Skip to content

Commit e9495e7

Browse files
authored
Merge pull request #4508 from Yancey1989/seqconcat_op
Add the sequence_concat operator.
2 parents 4b1f70d + d68122f commit e9495e7

File tree

4 files changed

+386
-0
lines changed

4 files changed

+386
-0
lines changed
Lines changed: 129 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,129 @@
1+
/* Copyright (c) 2016 PaddlePaddle Authors. All Rights Reserve.
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+
#include "paddle/operators/sequence_concat_op.h"
16+
17+
namespace paddle {
18+
namespace operators {
19+
20+
class SequenceConcatOp : public framework::OperatorWithKernel {
21+
public:
22+
using framework::OperatorWithKernel::OperatorWithKernel;
23+
24+
protected:
25+
void InferShape(framework::InferShapeContext* ctx) const override {
26+
PADDLE_ENFORCE(ctx->HasInputs("X"),
27+
"Inputs(X) of SequenceConcatOp should not be null.");
28+
PADDLE_ENFORCE(ctx->HasOutput("Out"),
29+
"Output(Out) of SequenceConcatOp should not be null.");
30+
const size_t level = static_cast<size_t>(ctx->Attrs().Get<int>("level"));
31+
const size_t axis = static_cast<size_t>(ctx->Attrs().Get<int>("axis"));
32+
PADDLE_ENFORCE(level == 0UL || level == 1UL,
33+
"The sequence_concat operator only accepts sequence "
34+
"or a nested sequence as its input.");
35+
auto ins_dims = ctx->GetInputsDim("X");
36+
framework::DDim out_dims = ins_dims[0];
37+
const size_t n = ins_dims.size();
38+
for (size_t i = 1; i < n; ++i) {
39+
out_dims[axis] += ins_dims[i][axis];
40+
}
41+
ctx->SetOutputDim("Out", out_dims);
42+
}
43+
};
44+
45+
class SequenceConcatOpMaker : public framework::OpProtoAndCheckerMaker {
46+
public:
47+
SequenceConcatOpMaker(framework::OpProto* proto,
48+
framework::OpAttrChecker* op_checker)
49+
: OpProtoAndCheckerMaker(proto, op_checker) {
50+
AddInput("X",
51+
"(A vector of LoDTensor), the input is a vector of LoDTensor, "
52+
"each of which is a variable-length sequence or nested sequence.")
53+
.AsDuplicable();
54+
AddOutput("Out",
55+
"(A LoDTensor), the variable-length output of "
56+
"sequence_concat Op.");
57+
AddAttr<int>("axis",
58+
"(int, default 0)"
59+
"The axis which the inputs will be joined with. "
60+
"If axis is 0, the inputs will be joined with LoD index.")
61+
.SetDefault(0);
62+
AddAttr<int>("level",
63+
"(int, default 0)"
64+
"The level at which the inputs will be joined. "
65+
"If the level is 0, the inputs will be joined at the nested "
66+
"sequence level. "
67+
"If the level is 1, the inputs will be joined at the "
68+
"sequence level. "
69+
"The level should be less than the level number of inputs.")
70+
.SetDefault(0);
71+
AddComment(R"DOC(
72+
The sequence_concat operator concatenates multiple LoDTensors.
73+
It only supports sequence (LoD Tensor with level number is 1)
74+
or a nested sequence (LoD tensor with level number is 2) as its input.
75+
- Case1:
76+
If the axis is other than 0(here, axis is 1 and level is 1),
77+
each input should have the same LoD information and the LoD
78+
information of the output keeps the same as the input.
79+
80+
LoD(x0) = {{0,2,4}, {0,1,2,3,4}}; Dims(x0) = (4,3,4)
81+
LoD(x1) = {{0,2,4}, {0,1,2,3,4}}; Dims(x1) = (4,4,4)
82+
LoD(Out) = {{0,2,4}, {0,1,2,3,4}}; Dims(Out) = (4,7,4)
83+
84+
- Case2:
85+
If the axis is 0(here, leve is 0), the inputs are concatenated along
86+
time steps, the LoD information of the output need to re-compute.
87+
88+
LoD(x0) = {{0,2,4}, {0,1,2,3,4}}; Dims(x0) = (4,3,4)
89+
LoD(x1) = {{0,3,5}, {0,1,2,3,5}}; Dims(x1) = (5,3,4)
90+
LoD(Out) = {{0,5,9}, {0,1,2,3,4,5,6,7,9}}; Dims(Out) = (9,3,4)
91+
92+
- Case3:
93+
If the axis is 0(here, level is 1).
94+
95+
LoD(x0) = {{0,2,4}, {0,1,2,3,4}}; Dims(x0) = (4,3,4)
96+
LoD(x1) = {{0,3,5}, {0,1,3,4,5}}; Dims(x1) = (5,3,4)
97+
LoD(Out) = {{0,5,9}, {0,2,5,7,9}}; Dims(Out) = (9,3,4)
98+
99+
NOTE: The levels of all the inputs should be the same.
100+
)DOC");
101+
}
102+
};
103+
104+
class SequenceConcatGradOp : public framework::OperatorWithKernel {
105+
public:
106+
using framework::OperatorWithKernel::OperatorWithKernel;
107+
108+
protected:
109+
void InferShape(framework::InferShapeContext* ctx) const override {
110+
PADDLE_ENFORCE(ctx->HasInput(framework::GradVarName("Out")),
111+
"The gradient of Out should not be null.");
112+
PADDLE_ENFORCE(ctx->HasOutputs(framework::GradVarName("X")),
113+
"The gradient of X should not be null.");
114+
ctx->SetOutputsDim(framework::GradVarName("X"), ctx->GetInputsDim("X"));
115+
}
116+
};
117+
118+
} // namespace operators
119+
} // namespace paddle
120+
121+
namespace ops = paddle::operators;
122+
REGISTER_OP(sequence_concat, ops::SequenceConcatOp, ops::SequenceConcatOpMaker,
123+
sequence_concat_grad, ops::SequenceConcatGradOp);
124+
REGISTER_OP_CPU_KERNEL(
125+
sequence_concat,
126+
ops::SequenceConcatOpKernel<paddle::platform::CPUPlace, float>);
127+
REGISTER_OP_CPU_KERNEL(
128+
sequence_concat_grad,
129+
ops::SequenceConcatGradOpKernel<paddle::platform::CPUPlace, float>);
Lines changed: 25 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,25 @@
1+
/* Copyright (c) 2016 PaddlePaddle Authors. All Rights Reserve.
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+
#define EIGEN_USE_GPU
16+
17+
#include "paddle/operators/sequence_concat_op.h"
18+
19+
namespace ops = paddle::operators;
20+
REGISTER_OP_GPU_KERNEL(
21+
sequence_concat,
22+
ops::SequenceConcatOpKernel<paddle::platform::GPUPlace, float>);
23+
REGISTER_OP_GPU_KERNEL(
24+
sequence_concat_grad,
25+
ops::SequenceConcatGradOpKernel<paddle::platform::GPUPlace, float>);
Lines changed: 155 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,155 @@
1+
/* Copyright (c) 2016 PaddlePaddle Authors. All Rights Reserve.
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+
#pragma once
16+
#include "paddle/framework/op_registry.h"
17+
#include "paddle/operators/strided_memcpy.h"
18+
19+
namespace paddle {
20+
namespace operators {
21+
22+
using Tensor = framework::Tensor;
23+
using LoDTensor = framework::LoDTensor;
24+
using LoD = framework::LoD;
25+
26+
template <typename T>
27+
LoD concatLoD(const std::vector<const T*> ins, const size_t axis,
28+
const size_t level) {
29+
auto out_lod = ins[0]->lod();
30+
const size_t n = ins.size();
31+
if (axis == 0UL) {
32+
for (size_t i = 1; i < n; ++i) {
33+
for (size_t j = 0; j < ins[i]->lod()[0].size(); ++j) {
34+
out_lod[0][j] += ins[i]->lod()[0][j];
35+
}
36+
37+
if (ins[0]->NumLevels() == 2) {
38+
for (size_t j = 1; j < ins[i]->lod()[1].size(); ++j) {
39+
if (level == 0UL) {
40+
out_lod[1].push_back(out_lod[1].back() + ins[i]->lod()[1][j] -
41+
ins[i]->lod()[1][j - 1]);
42+
} else if (level == 1UL) {
43+
out_lod[1][j] += ins[1]->lod()[1][j];
44+
}
45+
}
46+
}
47+
}
48+
}
49+
return out_lod;
50+
}
51+
52+
template <typename Place, typename T>
53+
class SequenceConcatOpKernel : public framework::OpKernel<T> {
54+
public:
55+
void Compute(const framework::ExecutionContext& ctx) const override {
56+
auto ins = ctx.MultiInput<LoDTensor>("X");
57+
auto* out = ctx.Output<LoDTensor>("Out");
58+
const size_t axis = static_cast<size_t>(ctx.Attr<int>("axis"));
59+
const size_t level = static_cast<size_t>(ctx.Attr<int>("level"));
60+
const size_t n = ins.size();
61+
62+
for (size_t i = 1; i < n; ++i) {
63+
PADDLE_ENFORCE_EQ(ins[0]->NumLevels(), ins[i]->NumLevels(),
64+
"The levels of all the input LoDTensors "
65+
"should be the same.");
66+
PADDLE_ENFORCE_EQ(ins[0]->dims().size(), ins[i]->dims().size(),
67+
"The dimension size of all the input LoDTensors "
68+
"should be the same.");
69+
70+
const size_t dims_size = ins[i]->dims().size();
71+
for (size_t j = 0; j < dims_size; ++j) {
72+
if (j == axis) continue;
73+
PADDLE_ENFORCE_EQ(ins[0]->dims()[j], ins[i]->dims()[j],
74+
"Except for the dimension of the specified "
75+
"axis along which all the inputs are concatenated, "
76+
"dimensions of all the other axises of the input "
77+
"LoDTensors should be the same.");
78+
}
79+
}
80+
PADDLE_ENFORCE_GT(ins[0]->NumLevels(), level,
81+
"The levels of all the input LoDTensors "
82+
"should be greater than the specify level");
83+
84+
out->mutable_data<T>(ctx.GetPlace());
85+
auto out_lod = concatLoD<LoDTensor>(ins, axis, level);
86+
out->set_lod(out_lod);
87+
88+
auto out_lod_level = out_lod[level];
89+
for (size_t i = 0; i < out_lod_level.size() - 1; ++i) {
90+
Tensor out_t = out->Slice<T>(static_cast<int>(out_lod_level[i]),
91+
static_cast<int>(out_lod_level[i + 1]));
92+
auto out_stride = framework::stride(out_t.dims());
93+
size_t offset = 0;
94+
95+
for (size_t j = 0; j < n; ++j) {
96+
auto in_lod_level = ins[j]->lod()[level];
97+
auto in_stride = framework::stride(ins[j]->dims());
98+
Tensor in_t = ins[j]->Slice<T>(static_cast<int>(in_lod_level[i]),
99+
static_cast<int>(in_lod_level[i + 1]));
100+
size_t axis_dim = in_t.dims()[axis];
101+
StridedMemcpy<T>(ctx.device_context(), in_t.data<T>(), in_stride,
102+
in_t.dims(), out_stride, out_t.data<T>() + offset);
103+
offset += axis_dim * in_stride[axis];
104+
}
105+
}
106+
}
107+
};
108+
109+
template <typename Place, typename T>
110+
class SequenceConcatGradOpKernel : public framework::OpKernel<T> {
111+
public:
112+
void Compute(const framework::ExecutionContext& ctx) const override {
113+
auto ins = ctx.MultiInput<framework::LoDTensor>("X");
114+
auto* out_grad =
115+
ctx.Input<framework::LoDTensor>(framework::GradVarName("Out"));
116+
auto x_grads =
117+
ctx.MultiOutput<framework::LoDTensor>(framework::GradVarName("X"));
118+
size_t axis = static_cast<size_t>(ctx.Attr<int>("axis"));
119+
size_t level = static_cast<size_t>(ctx.Attr<int>("level"));
120+
const size_t n = x_grads.size();
121+
122+
// Set Grad(X) LoD as X
123+
for (size_t i = 0; i < n; i++) {
124+
x_grads[i]->set_lod(ins[i]->lod());
125+
x_grads[i]->mutable_data<T>(ctx.GetPlace());
126+
}
127+
128+
auto out_lod = concatLoD<LoDTensor>(ins, axis, level);
129+
auto out_lod_level = out_lod[level];
130+
131+
for (size_t i = 0; i < out_lod_level.size() - 1; ++i) {
132+
Tensor out_grad_t =
133+
out_grad->Slice<T>(static_cast<int>(out_lod_level[i]),
134+
static_cast<int>(out_lod_level[i + 1]));
135+
auto out_grad_stride = framework::stride(out_grad_t.dims());
136+
size_t offset = 0;
137+
138+
for (size_t j = 0; j < n; ++j) {
139+
auto x_grad_lod_level = x_grads[j]->lod()[level];
140+
auto x_grad_stride = framework::stride(x_grads[j]->dims());
141+
Tensor x_grad_t =
142+
x_grads[j]->Slice<T>(static_cast<int>(x_grad_lod_level[i]),
143+
static_cast<int>(x_grad_lod_level[i + 1]));
144+
size_t axis_dim = x_grad_t.dims()[axis];
145+
StridedMemcpy<T>(ctx.device_context(), out_grad_t.data<T>() + offset,
146+
out_grad_stride, out_grad_t.dims(), x_grad_stride,
147+
x_grad_t.data<T>());
148+
offset += axis_dim * out_grad_stride[axis];
149+
}
150+
}
151+
}
152+
};
153+
154+
} // namespace operators
155+
} // namespace paddle
Lines changed: 77 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,77 @@
1+
import unittest
2+
import numpy as np
3+
from op_test import OpTest
4+
5+
6+
class TestConcatOp(OpTest):
7+
def set_data(self):
8+
# two level, batch size is 3
9+
x0 = np.random.random((4, 6, 3)).astype('float32')
10+
lod0 = [[0, 2, 4], [0, 1, 2, 3, 4]]
11+
x1 = np.random.random((4, 8, 3)).astype('float32')
12+
lod1 = [[0, 2, 4], [0, 1, 2, 3, 4]]
13+
axis = 1
14+
level = 1
15+
self.inputs = {'X': [('x0', (x0, lod0)), ('x1', (x1, lod1))]}
16+
self.attrs = {'axis': axis, 'level': level}
17+
outs = []
18+
for i in range(4):
19+
sub_x0 = x0[lod0[level][i]:lod0[level][i + 1], :]
20+
sub_x1 = x1[lod1[level][i]:lod1[level][i + 1], :]
21+
outs.append(np.concatenate((sub_x0, sub_x1), axis=axis))
22+
23+
self.outputs = {'Out': np.concatenate(outs, axis=0)}
24+
25+
def setUp(self):
26+
self.op_type = "sequence_concat"
27+
self.set_data()
28+
29+
def test_check_output(self):
30+
self.check_output()
31+
32+
def test_check_grad(self):
33+
self.check_grad(['x0'], 'Out')
34+
35+
36+
class TestConcatOpDiffLod(TestConcatOp):
37+
def set_data(self):
38+
# two level, batch size is 3
39+
x0 = np.random.random((4, 6, 3)).astype('float32')
40+
lod0 = [[0, 2, 4], [0, 1, 2, 3, 4]]
41+
x1 = np.random.random((5, 6, 3)).astype('float32')
42+
lod1 = [[0, 3, 5], [0, 1, 2, 3, 5]]
43+
axis = 0
44+
level = 1
45+
self.inputs = {'X': [('x0', (x0, lod0)), ('x1', (x1, lod1))]}
46+
self.attrs = {'axis': axis, 'level': level}
47+
outs = []
48+
for i in range(4):
49+
sub_x0 = x0[lod0[level][i]:lod0[level][i + 1], :]
50+
sub_x1 = x1[lod1[level][i]:lod1[level][i + 1], :]
51+
outs.append(np.concatenate((sub_x0, sub_x1), axis=axis))
52+
53+
self.outputs = {'Out': np.concatenate(outs, axis=0)}
54+
55+
56+
class TestConcatOpLevelZero(TestConcatOp):
57+
def set_data(self):
58+
# two level, batch size is 3
59+
x0 = np.random.random((4, 3, 4)).astype('float32')
60+
lod0 = [[0, 2, 4], [0, 1, 2, 3, 4]]
61+
x1 = np.random.random((5, 3, 4)).astype('float32')
62+
lod1 = [[0, 3, 5], [0, 1, 3, 4, 5]]
63+
axis = 0
64+
level = 0
65+
self.inputs = {'X': [('x0', (x0, lod0)), ('x1', (x1, lod1))]}
66+
self.attrs = {'axis': axis, 'level': level}
67+
outs = []
68+
for i in range(2):
69+
sub_x0 = x0[lod0[level][i]:lod0[level][i + 1], :]
70+
sub_x1 = x1[lod1[level][i]:lod1[level][i + 1], :]
71+
outs.append(np.concatenate((sub_x0, sub_x1), axis=axis))
72+
73+
self.outputs = {'Out': np.concatenate(outs, axis=0)}
74+
75+
76+
if __name__ == '__main__':
77+
unittest.main()

0 commit comments

Comments
 (0)