-
Notifications
You must be signed in to change notification settings - Fork 487
/
tabulate_multi_device.cc
263 lines (246 loc) · 11.5 KB
/
tabulate_multi_device.cc
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
#include "custom_op.h"
#include "tabulate.h"
REGISTER_OP("TabulateFusion")
.Attr("T: {float, double} = DT_DOUBLE")
.Input("table: T")
.Input("table_info: T")
.Input("em_x: T")
.Input("em: T")
.Attr("last_layer_size: int")
.Output("descriptor: T");
REGISTER_OP("TabulateFusionGrad")
.Attr("T: {float, double} = DT_DOUBLE")
.Input("table: T")
.Input("table_info: T")
.Input("em_x: T")
.Input("em: T")
.Input("dy: T")
.Input("descriptor: T")
.Output("dy_dem_x: T")
.Output("dy_dem: T");
REGISTER_OP("TabulateFusionGradGrad")
.Attr("T: {float, double}")
.Input("table: T")
.Input("table_info: T")
.Input("em_x: T")
.Input("em: T")
.Input("dz_dy_dem_x: T")
.Input("dz_dy_dem: T")
.Input("descriptor: T")
.Output("dz_dy: T");
template<typename Device, typename FPTYPE>
class TabulateFusionOp : public OpKernel {
public:
explicit TabulateFusionOp(OpKernelConstruction* context) : OpKernel(context) {
OP_REQUIRES_OK(context, context->GetAttr("last_layer_size", &last_layer_size));
}
void Compute(OpKernelContext* context) override {
deepmd::safe_compute(context, [this](OpKernelContext* context) {this->_Compute(context);});
}
void _Compute(OpKernelContext* context) {
// Grab the input tensor
int context_input_index = 0;
const Tensor& table_tensor = context->input(context_input_index++);
const Tensor& table_info_tensor = context->input(context_input_index++);
const Tensor& em_x_tensor = context->input(context_input_index++);
const Tensor& em_tensor = context->input(context_input_index++);
// set size of the sample
OP_REQUIRES (context, (table_tensor.shape().dims() == 2), errors::InvalidArgument ("Dim of table should be 2"));
OP_REQUIRES (context, (em_x_tensor.shape().dims() == 2), errors::InvalidArgument ("Dim of input should be 2"));
OP_REQUIRES (context, (em_tensor.shape().dims() == 3), errors::InvalidArgument ("Dim of input should be 3"));
TensorShape descriptor_shape;
descriptor_shape.AddDim (em_tensor.shape().dim_size(0));
descriptor_shape.AddDim (4); // TODO: be careful here;
descriptor_shape.AddDim (last_layer_size);
int context_output_index = 0;
Tensor* descriptor_tensor = NULL;
OP_REQUIRES_OK(context, context->allocate_output(
context_output_index++,
descriptor_shape,
&descriptor_tensor));
DeviceFunctor() (
device,
context->eigen_device<Device>()
);
// flat the tensors
FPTYPE * descriptor = descriptor_tensor->flat<FPTYPE>().data();
const FPTYPE * table = table_tensor.flat<FPTYPE>().data();
const FPTYPE * table_info = table_info_tensor.flat<FPTYPE>().data();
const FPTYPE * em_x = em_x_tensor.flat<FPTYPE>().data();
const FPTYPE * em = em_tensor.flat<FPTYPE>().data();
const int nloc = em_tensor.shape().dim_size(0);
const int nnei = em_tensor.shape().dim_size(1);
if (device == "GPU") {
#if GOOGLE_CUDA
deepmd::tabulate_fusion_gpu_cuda(
descriptor,
table, table_info, em_x, em, nloc, nnei, last_layer_size);
#endif // GOOGLE_CUDA
#if TENSORFLOW_USE_ROCM
deepmd::tabulate_fusion_gpu_rocm(
descriptor,
table, table_info, em_x, em, nloc, nnei, last_layer_size);
#endif // TENSORFLOW_USE_ROCM
}
else if (device == "CPU") {
deepmd::tabulate_fusion_cpu(
descriptor,
table, table_info, em_x, em, nloc, nnei, last_layer_size);
}
}
private:
int last_layer_size;
std::string device;
};
template<typename Device, typename FPTYPE>
class TabulateFusionGradOp : public OpKernel {
public:
explicit TabulateFusionGradOp(OpKernelConstruction* context) : OpKernel(context) {}
void Compute(OpKernelContext* context) override {
deepmd::safe_compute(context, [this](OpKernelContext* context) {this->_Compute(context);});
}
void _Compute(OpKernelContext* context) {
// Grab the input tensor
int context_input_index = 0;
const Tensor& table_tensor = context->input(context_input_index++);
const Tensor& table_info_tensor = context->input(context_input_index++);
const Tensor& em_x_tensor = context->input(context_input_index++);
const Tensor& em_tensor = context->input(context_input_index++);
const Tensor& dy_tensor = context->input(context_input_index++);
const Tensor& descriptor_tensor = context->input(context_input_index++);
// set size of the sample
OP_REQUIRES (context, (dy_tensor.shape().dims() == 3), errors::InvalidArgument ("Dim of table should be 3"));
int context_output_index = 0;
Tensor* dy_dem_x_tensor = NULL;
OP_REQUIRES_OK(context, context->allocate_output(
context_output_index++,
em_x_tensor.shape(),
&dy_dem_x_tensor));
Tensor* dy_dem_tensor = NULL;
OP_REQUIRES_OK(context, context->allocate_output(
context_output_index++,
em_tensor.shape(),
&dy_dem_tensor));
DeviceFunctor() (
device,
context->eigen_device<Device>()
);
// flat the tensors
FPTYPE * dy_dem_x = dy_dem_x_tensor->flat<FPTYPE>().data();
FPTYPE * dy_dem = dy_dem_tensor->flat<FPTYPE>().data();
const FPTYPE * descriptor = descriptor_tensor.flat<FPTYPE>().data();
const FPTYPE * table = table_tensor.flat<FPTYPE>().data();
const FPTYPE * table_info = table_info_tensor.flat<FPTYPE>().data();
const FPTYPE * em_x = em_x_tensor.flat<FPTYPE>().data();
const FPTYPE * em = em_tensor.flat<FPTYPE>().data();
const FPTYPE * dy = dy_tensor.flat<FPTYPE>().data();
const int nloc = em_tensor.shape().dim_size(0);
const int nnei = em_tensor.shape().dim_size(1);
const int last_layer_size = descriptor_tensor.shape().dim_size(2);
if (device == "GPU") {
#if GOOGLE_CUDA
deepmd::tabulate_fusion_grad_gpu_cuda(
dy_dem_x, dy_dem,
table, table_info, em_x, em, dy, nloc, nnei, last_layer_size);
#endif // GOOGLE_CUDA
#if TENSORFLOW_USE_ROCM
deepmd::tabulate_fusion_grad_gpu_rocm(
dy_dem_x, dy_dem,
table, table_info, em_x, em, dy, nloc, nnei, last_layer_size);
#endif // TENSORFLOW_USE_ROCM
}
else if (device == "CPU") {
deepmd::tabulate_fusion_grad_cpu(
dy_dem_x, dy_dem,
table, table_info, em_x, em, dy, nloc, nnei, last_layer_size);
}
}
private:
std::string device;
};
template<typename Device, typename FPTYPE>
class TabulateFusionGradGradOp : public OpKernel {
public:
explicit TabulateFusionGradGradOp(OpKernelConstruction* context) : OpKernel(context) {}
void Compute(OpKernelContext* context) override {
// Grab the input tensor
int context_input_index = 0;
const Tensor& table_tensor = context->input(context_input_index++);
const Tensor& table_info_tensor = context->input(context_input_index++);
const Tensor& em_x_tensor = context->input(context_input_index++);
const Tensor& em_tensor = context->input(context_input_index++);
const Tensor& dz_dy_dem_x_tensor = context->input(context_input_index++);
const Tensor& dz_dy_dem_tensor = context->input(context_input_index++);
const Tensor& descriptor_tensor = context->input(context_input_index++);
// set size of the sample
OP_REQUIRES (context, (dz_dy_dem_x_tensor.shape().dims() == 2), errors::InvalidArgument ("Dim of input should be 2"));
OP_REQUIRES (context, (dz_dy_dem_tensor.shape().dims() == 3), errors::InvalidArgument ("Dim of input should be 3"));
int context_output_index = 0;
Tensor* dz_dy_tensor = NULL;
OP_REQUIRES_OK(context, context->allocate_output(
context_output_index++,
descriptor_tensor.shape(),
&dz_dy_tensor));
DeviceFunctor() (
device,
context->eigen_device<Device>()
);
// flat the tensors
FPTYPE * dz_dy = dz_dy_tensor->flat<FPTYPE>().data();
const FPTYPE * table = table_tensor.flat<FPTYPE>().data();
const FPTYPE * table_info = table_info_tensor.flat<FPTYPE>().data();
const FPTYPE * em_x = em_x_tensor.flat<FPTYPE>().data();
const FPTYPE * em = em_tensor.flat<FPTYPE>().data();
const FPTYPE * dz_dy_dem_x = dz_dy_dem_x_tensor.flat<FPTYPE>().data();
const FPTYPE * dz_dy_dem = dz_dy_dem_tensor.flat<FPTYPE>().data();
const int nloc = em_tensor.shape().dim_size(0);
const int nnei = em_tensor.shape().dim_size(1);
const int last_layer_size = descriptor_tensor.shape().dim_size(2);
if (device == "GPU") {
#if GOOGLE_CUDA
deepmd::tabulate_fusion_grad_grad_gpu_cuda(
dz_dy,
table, table_info, em_x, em, dz_dy_dem_x, dz_dy_dem, nloc, nnei, last_layer_size);
#endif // GOOGLE_CUDA
#if TENSORFLOW_USE_ROCM
deepmd::tabulate_fusion_grad_grad_gpu_rocm(
dz_dy,
table, table_info, em_x, em, dz_dy_dem_x, dz_dy_dem, nloc, nnei, last_layer_size);
#endif // TENSORFLOW_USE_ROCM
OP_REQUIRES (context, (last_layer_size <= 1024), errors::InvalidArgument ("In the process of model compression, the size of the last layer of embedding net must be less than 1024!"));
}
else if (device == "CPU") {
deepmd::tabulate_fusion_grad_grad_cpu(
dz_dy,
table, table_info, em_x, em, dz_dy_dem_x, dz_dy_dem, nloc, nnei, last_layer_size);
}
}
private:
std::string device;
};
#define REGISTER_CPU(T) \
REGISTER_KERNEL_BUILDER( \
Name("TabulateFusion").Device(DEVICE_CPU).TypeConstraint<T>("T").HostMemory("table_info"), \
TabulateFusionOp<CPUDevice, T>); \
REGISTER_KERNEL_BUILDER( \
Name("TabulateFusionGrad").Device(DEVICE_CPU).TypeConstraint<T>("T").HostMemory("table_info"), \
TabulateFusionGradOp<CPUDevice, T>); \
REGISTER_KERNEL_BUILDER( \
Name("TabulateFusionGradGrad").Device(DEVICE_CPU).TypeConstraint<T>("T").HostMemory("table_info"), \
TabulateFusionGradGradOp<CPUDevice, T>);
REGISTER_CPU(float);
REGISTER_CPU(double);
#if GOOGLE_CUDA || TENSORFLOW_USE_ROCM
#define REGISTER_GPU(T) \
REGISTER_KERNEL_BUILDER( \
Name("TabulateFusion").Device(DEVICE_GPU).TypeConstraint<T>("T").HostMemory("table_info"), \
TabulateFusionOp<GPUDevice, T>); \
REGISTER_KERNEL_BUILDER( \
Name("TabulateFusionGrad").Device(DEVICE_GPU).TypeConstraint<T>("T").HostMemory("table_info"), \
TabulateFusionGradOp<GPUDevice, T>); \
REGISTER_KERNEL_BUILDER( \
Name("TabulateFusionGradGrad").Device(DEVICE_GPU).TypeConstraint<T>("T").HostMemory("table_info"), \
TabulateFusionGradGradOp<GPUDevice, T>);
REGISTER_GPU(float);
REGISTER_GPU(double);
#endif // GOOGLE_CUDA || TENSORFLOW_USE_ROCM