• Home
  • Line#
  • Scopes#
  • Navigate#
  • Raw
  • Download
1 /* Copyright 2021 The TensorFlow Authors. All Rights Reserved.
2 
3 Licensed under the Apache License, Version 2.0 (the "License");
4 you may not use this file except in compliance with the License.
5 You may obtain a copy of the License at
6 
7     http://www.apache.org/licenses/LICENSE-2.0
8 
9 Unless required by applicable law or agreed to in writing, software
10 distributed under the License is distributed on an "AS IS" BASIS,
11 WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
12 See the License for the specific language governing permissions and
13 limitations under the License.
14 ==============================================================================*/
15 
16 #include "tensorflow/compiler/xla/service/gpu/nvptx_compiler.h"
17 
18 #include "tensorflow/compiler/xla/service/buffer_assignment.h"
19 #include "tensorflow/compiler/xla/service/hlo_instruction.h"
20 #include "tensorflow/compiler/xla/service/hlo_parser.h"
21 #include "tensorflow/compiler/xla/status_macros.h"
22 #include "tensorflow/compiler/xla/tests/hlo_test_base.h"
23 #include "tensorflow/compiler/xla/util.h"
24 
25 namespace xla {
26 namespace gpu {
27 
28 using NVPTXCompilerTest = HloTestBase;
29 
TEST_F(NVPTXCompilerTest,AllReducePerformedInplace)30 TEST_F(NVPTXCompilerTest, AllReducePerformedInplace) {
31   const absl::string_view hlo_string = R"(
32 HloModule Module, input_output_alias={ {}: (0, {}, may-alias) }
33 
34 summit {
35   lhs = f32[] parameter(0)
36   rhs = f32[] parameter(1)
37   ROOT add = f32[] add(lhs, rhs)
38 }
39 
40 ENTRY entry {
41   param0 = f32[128] parameter(0)
42   ROOT allreduce = f32[128] all-reduce(param0),
43     replica_groups={}, to_apply=summit
44 }
45 )";
46   TF_ASSERT_OK_AND_ASSIGN(std::unique_ptr<HloModule> module,
47                           ParseAndReturnVerifiedModule(hlo_string));
48 
49   NVPTXCompiler compiler;
50   TF_ASSERT_OK_AND_ASSIGN(auto buffer_assignment,
51                           compiler.AssignBuffers(module.get()));
52 
53   HloInstruction* all_reduce = module->entry_computation()->root_instruction();
54   EXPECT_TRUE(buffer_assignment->SharesTopLevelSlice(all_reduce,
55                                                      all_reduce->operand(0)));
56 }
57 
TEST_F(NVPTXCompilerTest,AllReducePerformedInplaceTwoOperands)58 TEST_F(NVPTXCompilerTest, AllReducePerformedInplaceTwoOperands) {
59   const absl::string_view hlo_string = R"(
60 HloModule Module,
61   input_output_alias={ {0}: (0, {}, may-alias), {1}: (1, {}, may-alias) }
62 
63 summit {
64   lhs = f32[] parameter(0)
65   rhs = f32[] parameter(1)
66   ROOT add = f32[] add(lhs, rhs)
67 }
68 
69 ENTRY entry {
70   param0 = f32[128] parameter(0)
71   param1 = f32[128] parameter(1)
72   ROOT allreduce = (f32[128], f32[128]) all-reduce(param0, param1),
73     replica_groups={}, to_apply=summit
74 }
75 )";
76   TF_ASSERT_OK_AND_ASSIGN(std::unique_ptr<HloModule> module,
77                           ParseAndReturnVerifiedModule(hlo_string));
78 
79   NVPTXCompiler compiler;
80   TF_ASSERT_OK_AND_ASSIGN(auto buffer_assignment,
81                           compiler.AssignBuffers(module.get()));
82 
83   HloInstruction* all_reduce = module->entry_computation()->root_instruction();
84   EXPECT_TRUE(buffer_assignment->SharesSliceAtIndex(
85       all_reduce, {0}, all_reduce->operand(0), {}));
86   EXPECT_TRUE(buffer_assignment->SharesSliceAtIndex(
87       all_reduce, {1}, all_reduce->operand(1), {}));
88 }
89 
90 }  // namespace gpu
91 }  // namespace xla
92