Skip to content

Commit f276006

Browse files
authored
Merge pull request #12694 from JiayiFeng/dev_op_tensor_support
Make cross_entropy_op supporting tensor
2 parents a197737 + d6b5302 commit f276006

File tree

6 files changed

+184
-58
lines changed

6 files changed

+184
-58
lines changed

paddle/fluid/framework/tensor.cc

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -112,5 +112,6 @@ Tensor& Tensor::Resize(const DDim& dims) {
112112
const DDim& Tensor::dims() const { return dims_; }
113113

114114
int64_t Tensor::numel() const { return product(dims_); }
115+
115116
} // namespace framework
116117
} // namespace paddle

paddle/fluid/framework/tensor_impl.h

Lines changed: 8 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -59,6 +59,14 @@ inline T* Tensor::mutable_data(platform::Place place) {
5959
}
6060

6161
inline Tensor ReshapeToMatrix(const Tensor& src, int num_col_dims) {
62+
int rank = src.dims().size();
63+
PADDLE_ENFORCE_GE(
64+
rank, 2,
65+
"'ReshapeToMatrix()' is only used for flatten high rank "
66+
"tensors to matrixs. Can not be used in reshaping vectors.");
67+
if (rank == 2) {
68+
return src;
69+
}
6270
Tensor res;
6371
res.ShareDataWith(src);
6472
res.Resize(flatten_to_2d(src.dims(), num_col_dims));

paddle/fluid/operators/cross_entropy_op.cc

Lines changed: 52 additions & 37 deletions
Original file line numberDiff line numberDiff line change
@@ -28,23 +28,26 @@ class CrossEntropyOp : public framework::OperatorWithKernel {
2828

2929
auto x_dims = ctx->GetInputDim("X");
3030
auto label_dims = ctx->GetInputDim("Label");
31-
PADDLE_ENFORCE_EQ(x_dims.size(), 2UL, "Input(X)'s rank should be 2.");
32-
PADDLE_ENFORCE_EQ(label_dims.size(), 2UL,
33-
"Input(Label)'s rank should be 2.");
34-
PADDLE_ENFORCE_EQ(x_dims[0], label_dims[0],
35-
"The 1st dimension of Input(X) and Input(Label) should "
36-
"be equal.");
31+
int rank = x_dims.size();
32+
PADDLE_ENFORCE_EQ(rank, label_dims.size(),
33+
"Input(X) and Input(Label) shall have the same rank.");
34+
PADDLE_ENFORCE_EQ(framework::slice_ddim(x_dims, 0, rank - 1),
35+
framework::slice_ddim(label_dims, 0, rank - 1),
36+
"Input(X) and Input(Label) shall have the same shape "
37+
"except the last dimension.");
3738
if (ctx->Attrs().Get<bool>("soft_label")) {
38-
PADDLE_ENFORCE_EQ(x_dims[1], label_dims[1],
39-
"If Attr(soft_label) == true, the 2nd dimension of "
39+
PADDLE_ENFORCE_EQ(x_dims[rank - 1], label_dims[rank - 1],
40+
"If Attr(soft_label) == true, the last dimension of "
4041
"Input(X) and Input(Label) should be equal.");
4142
} else {
42-
PADDLE_ENFORCE_EQ(label_dims[1], 1UL,
43-
"If Attr(softLabel) == false, the 2nd dimension of "
43+
PADDLE_ENFORCE_EQ(label_dims[rank - 1], 1UL,
44+
"If Attr(softLabel) == false, the last dimension of "
4445
"Input(Label) should be 1.");
4546
}
4647

47-
ctx->SetOutputDim("Y", {x_dims[0], 1});
48+
auto y_dims = x_dims;
49+
y_dims[rank - 1] = 1;
50+
ctx->SetOutputDim("Y", y_dims);
4851
ctx->ShareLoD("X", /*->*/ "Y");
4952
}
5053

@@ -74,24 +77,28 @@ class CrossEntropyGradientOp : public framework::OperatorWithKernel {
7477
auto x_dims = ctx->GetInputDim("X");
7578
auto label_dims = ctx->GetInputDim("Label");
7679
auto dy_dims = ctx->GetInputDim(framework::GradVarName("Y"));
77-
PADDLE_ENFORCE_EQ(x_dims.size(), 2, "Input(X)'s rank should be 2.");
78-
PADDLE_ENFORCE_EQ(dy_dims.size(), 2, "Input(Y@Grad)'s rank should be 2.");
79-
PADDLE_ENFORCE_EQ(label_dims.size(), 2, "Input(Label)'s rank should be 2.");
80-
PADDLE_ENFORCE_EQ(x_dims[0], label_dims[0],
81-
"The 1st dimension of Input(X) and Input(Label) should "
82-
"be equal.");
83-
PADDLE_ENFORCE_EQ(x_dims[0], dy_dims[0],
84-
"The 1st dimension of Input(X) and Input(Y@Grad) should "
85-
"be equal.");
86-
PADDLE_ENFORCE_EQ(dy_dims[1], 1,
87-
"The 2nd dimension of Input(Y@Grad) should be 1.");
80+
int rank = x_dims.size();
81+
PADDLE_ENFORCE_EQ(dy_dims.size(), rank,
82+
"Input(Y@Grad) and Input(X) should have the same rank.");
83+
PADDLE_ENFORCE_EQ(label_dims.size(), rank,
84+
"Input(Label) and Input(X) should have the same rank.");
85+
PADDLE_ENFORCE_EQ(framework::slice_ddim(x_dims, 0, rank - 1),
86+
framework::slice_ddim(label_dims, 0, rank - 1),
87+
"The Input(X) and Input(Label) should have the same "
88+
"shape except the last dimension.");
89+
PADDLE_ENFORCE_EQ(framework::slice_ddim(x_dims, 0, rank - 1),
90+
framework::slice_ddim(dy_dims, 0, rank - 1),
91+
"The Input(X) and Input(Y@Grad) should have the same "
92+
"shape except the last dimension.");
93+
PADDLE_ENFORCE_EQ(dy_dims[rank - 1], 1,
94+
"The last dimension of Input(Y@Grad) should be 1.");
8895
if (ctx->Attrs().Get<bool>("soft_label")) {
89-
PADDLE_ENFORCE_EQ(x_dims[1], label_dims[1],
90-
"When Attr(soft_label) == true, the 2nd dimension of "
96+
PADDLE_ENFORCE_EQ(x_dims[rank - 1], label_dims[rank - 1],
97+
"When Attr(soft_label) == true, the last dimension of "
9198
"Input(X) and Input(Label) should be equal.");
9299
} else {
93-
PADDLE_ENFORCE_EQ(label_dims[1], 1,
94-
"When Attr(soft_label) == false, the 2nd dimension of "
100+
PADDLE_ENFORCE_EQ(label_dims[rank - 1], 1,
101+
"When Attr(soft_label) == false, the last dimension of "
95102
"Input(Label) should be 1.");
96103
}
97104
ctx->SetOutputDim(framework::GradVarName("X"), x_dims);
@@ -113,25 +120,33 @@ class CrossEntropyOpMaker : public framework::OpProtoAndCheckerMaker {
113120
public:
114121
void Make() override {
115122
AddInput("X",
116-
"(Tensor, default Tensor<float>), a 2-D tensor with shape [N x D],"
117-
" where N is the batch size and D is the number of classes. "
118-
"This input is a probability computed by the previous operator, "
119-
"which is almost always the result of a softmax operator.");
120-
AddInput("Label",
121-
"(Tensor), the ground truth which is a 2-D tensor. When "
122-
"soft_label is set to false, Label is a Tensor<int64> with shape "
123-
"[N x 1]. When soft_label is set to true, Label is a "
124-
"Tensor<float/double> with shape [N x D].");
123+
"(Tensor, default Tensor<float>), a tensor whose last dimension "
124+
"size is equal to the number of classes. This input is a "
125+
"probability computed by the previous operator, which is almost "
126+
"always the result of a softmax operator.");
127+
AddInput(
128+
"Label",
129+
"(Tensor), the tensor which represents the ground truth. It has the "
130+
"same shape with 'X' except the last dimension. When soft_label is set "
131+
"to false, the last dimension size is 1; when soft_label is set to "
132+
"true, the last dimension size is equal to the number of classes.");
125133
AddOutput("Y",
126-
"(Tensor, default Tensor<float>), a 2-D tensor with shape "
127-
"[N x 1]. The cross entropy loss.");
134+
"(Tensor, default Tensor<float>), a tensor whose shape is same "
135+
"with 'X' except that the last dimension size is 1. It "
136+
"represents the cross entropy loss.");
128137
AddAttr<bool>("soft_label",
129138
"(bool, default false), a flag indicating whether to "
130139
"interpretate the given labels as soft labels.")
131140
.SetDefault(false);
132141
AddComment(R"DOC(
133142
CrossEntropy Operator.
134143
144+
The input 'X' and 'Label' will first be logically flattened to 2-D matrixs.
145+
The matrix's second dimension(row length) is as same as the original last
146+
dimension, and the first dimension(column length) is the product of all other
147+
original dimensions. Then the softmax computation will take palce on each raw
148+
of flattened matrixs.
149+
135150
It supports both standard cross-entropy and soft-label cross-entropy loss
136151
computation.
137152
1) One-hot cross-entropy:

paddle/fluid/operators/cross_entropy_op.h

Lines changed: 11 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -33,8 +33,13 @@ class CrossEntropyOpKernel : public framework::OpKernel<T> {
3333
auto* y = ctx.Output<Tensor>("Y");
3434
y->mutable_data<T>(ctx.GetPlace());
3535

36+
int rank = x->dims().size();
37+
Tensor x_2d = framework::ReshapeToMatrix(*x, rank - 1);
38+
Tensor labels_2d = framework::ReshapeToMatrix(*labels, rank - 1);
39+
Tensor y_2d = framework::ReshapeToMatrix(*y, rank - 1);
40+
3641
math::CrossEntropyFunctor<DeviceContext, T>()(
37-
ctx.template device_context<DeviceContext>(), y, x, labels,
42+
ctx.template device_context<DeviceContext>(), &y_2d, &x_2d, &labels_2d,
3843
ctx.Attr<bool>("soft_label"));
3944
}
4045
};
@@ -98,9 +103,12 @@ class CrossEntropyGradientOpKernel : public framework::OpKernel<T> {
98103
auto* dy = ctx.Input<Tensor>(framework::GradVarName("Y"));
99104
auto* label = ctx.Input<Tensor>("Label");
100105
auto* dx = ctx.Output<Tensor>(framework::GradVarName("X"));
101-
auto* dx_data = dx->mutable_data<T>(ctx.GetPlace());
106+
T* dx_data = dx->mutable_data<T>(ctx.GetPlace());
102107

103-
int64_t class_num = x->dims()[1];
108+
// Following computation only depends on the last dimension size. So it's
109+
// unnecessary to convert tensors to 2-D views.
110+
int rank = x->dims().size();
111+
int64_t class_num = x->dims()[rank - 1];
104112
if (ctx.Attr<bool>("soft_label")) {
105113
XeSoftlabelGradFunctor<T> functor(dx_data, dy->data<T>(), x->data<T>(),
106114
label->data<T>(),

paddle/fluid/operators/softmax_op.h

Lines changed: 10 additions & 18 deletions
Original file line numberDiff line numberDiff line change
@@ -31,16 +31,12 @@ class SoftmaxKernel : public framework::OpKernel<T> {
3131
// allocate memory on device.
3232
Out->mutable_data<T>(context.GetPlace());
3333

34-
auto dims = X->dims();
35-
auto flattened_dims = framework::flatten_to_2d(dims, dims.size() - 1);
36-
framework::LoDTensor flattened_x;
37-
framework::LoDTensor flattened_out;
38-
flattened_x.ShareDataWith(*X).Resize(flattened_dims);
39-
flattened_out.ShareDataWith(*Out).Resize(flattened_dims);
34+
int rank = X->dims().size();
35+
Tensor X_2d = framework::ReshapeToMatrix(*X, rank - 1);
36+
Tensor Out_2d = framework::ReshapeToMatrix(*Out, rank - 1);
4037

4138
math::SoftmaxFunctor<DeviceContext, T>()(
42-
context.template device_context<DeviceContext>(), &flattened_x,
43-
&flattened_out);
39+
context.template device_context<DeviceContext>(), &X_2d, &Out_2d);
4440
}
4541
};
4642

@@ -55,18 +51,14 @@ class SoftmaxGradKernel : public framework::OpKernel<T> {
5551
// allocate memory on device.
5652
dX->mutable_data<T>(context.GetPlace());
5753

58-
auto dims = Out->dims();
59-
auto flattened_dims = framework::flatten_to_2d(dims, dims.size() - 1);
60-
framework::LoDTensor flattened_out;
61-
framework::LoDTensor flattened_d_out;
62-
framework::LoDTensor flattened_d_x;
63-
flattened_out.ShareDataWith(*Out).Resize(flattened_dims);
64-
flattened_d_out.ShareDataWith(*dOut).Resize(flattened_dims);
65-
flattened_d_x.ShareDataWith(*dX).Resize(flattened_dims);
54+
int rank = Out->dims().size();
55+
Tensor Out_2d = framework::ReshapeToMatrix(*Out, rank - 1);
56+
Tensor dOut_2d = framework::ReshapeToMatrix(*dOut, rank - 1);
57+
Tensor dX_2d = framework::ReshapeToMatrix(*dX, rank - 1);
6658

6759
math::SoftmaxGradFunctor<DeviceContext, T>()(
68-
context.template device_context<DeviceContext>(), &flattened_out,
69-
&flattened_d_out, &flattened_d_x);
60+
context.template device_context<DeviceContext>(), &Out_2d, &dOut_2d,
61+
&dX_2d);
7062
}
7163
};
7264

python/paddle/fluid/tests/unittests/test_cross_entropy_op.py

Lines changed: 102 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -105,5 +105,107 @@ def test_check_grad(self):
105105
["X"], "Y", max_relative_error=0.05, numeric_grad_delta=0.001)
106106

107107

108+
class TestCrossEntropyOp4(OpTest):
109+
"""Test high rank tensor cross-entropy with discrete one-hot labels.
110+
"""
111+
112+
def setUp(self):
113+
self.op_type = "cross_entropy"
114+
shape = [10, 2, 4]
115+
ins_num = np.prod(np.array(shape))
116+
class_num = 10
117+
118+
X_2d = randomize_probability(ins_num, class_num, dtype='float64')
119+
120+
label_2d = np.random.randint(0, class_num, (ins_num, 1), dtype="int64")
121+
cross_entropy_2d = np.asmatrix(
122+
[[-np.log(X_2d[i][label_2d[i][0]])] for i in range(X_2d.shape[0])],
123+
dtype="float64")
124+
125+
X = X_2d.reshape(shape + [class_num])
126+
label = label_2d.reshape(shape + [1])
127+
cross_entropy = np.array(cross_entropy_2d).reshape(shape + [1])
128+
129+
self.inputs = {"X": X, "Label": label}
130+
self.outputs = {"Y": cross_entropy}
131+
self.attrs = {"soft_label": False}
132+
133+
def test_check_output(self):
134+
self.check_output()
135+
136+
def test_check_grad(self):
137+
self.check_grad(["X"], "Y", numeric_grad_delta=0.001)
138+
139+
140+
class TestCrossEntropyOp5(OpTest):
141+
"""Test high rank tensor cross-entropy with vectorized soft labels.
142+
"""
143+
144+
def setUp(self):
145+
self.op_type = "cross_entropy"
146+
shape = [4, 3]
147+
ins_num = np.prod(np.array(shape))
148+
class_num = 37
149+
150+
X_2d = randomize_probability(ins_num, class_num)
151+
label_2d = np.random.uniform(0.1, 1.0,
152+
[ins_num, class_num]).astype("float32")
153+
label_2d /= label_2d.sum(axis=1, keepdims=True)
154+
cross_entropy_2d = (-label_2d * np.log(X_2d)).sum(
155+
axis=1, keepdims=True).astype("float32")
156+
157+
X = X_2d.reshape(shape + [class_num])
158+
label = label_2d.reshape(shape + [class_num])
159+
cross_entropy = np.array(cross_entropy_2d).reshape(shape + [1])
160+
161+
self.inputs = {"X": X, "Label": label}
162+
self.outputs = {"Y": cross_entropy}
163+
self.attrs = {"soft_label": True}
164+
165+
def test_check_output(self):
166+
self.check_output()
167+
168+
def test_check_grad(self):
169+
self.check_grad(
170+
["X"], "Y", max_relative_error=0.05, numeric_grad_delta=0.001)
171+
172+
173+
class TestCrossEntropyOp6(OpTest):
174+
"""Test high rank tensor cross-entropy with vectorized one-hot representation of labels.
175+
"""
176+
177+
def setUp(self):
178+
self.op_type = "cross_entropy"
179+
shape = [4, 3, 2]
180+
ins_num = np.prod(np.array(shape))
181+
class_num = 17
182+
183+
X_2d = randomize_probability(ins_num, class_num)
184+
label_index_2d = np.random.randint(
185+
0, class_num, (ins_num), dtype="int32")
186+
label_2d = np.zeros(X_2d.shape)
187+
label_2d[np.arange(ins_num), label_index_2d] = 1
188+
189+
cross_entropy_2d = np.asmatrix(
190+
[[-np.log(X_2d[i][label_index_2d[i]])]
191+
for i in range(X_2d.shape[0])],
192+
dtype="float32")
193+
194+
X = X_2d.reshape(shape + [class_num])
195+
label = label_2d.reshape(shape + [class_num])
196+
cross_entropy = np.array(cross_entropy_2d).reshape(shape + [1])
197+
198+
self.inputs = {"X": X, "Label": label.astype(np.float32)}
199+
self.outputs = {"Y": cross_entropy}
200+
self.attrs = {"soft_label": True}
201+
202+
def test_check_output(self):
203+
self.check_output()
204+
205+
def test_check_grad(self):
206+
self.check_grad(
207+
["X"], "Y", max_relative_error=0.05, numeric_grad_delta=0.001)
208+
209+
108210
if __name__ == "__main__":
109211
unittest.main()

0 commit comments

Comments
 (0)