Skip to content

Commit d43932c

Browse files
authored
Merge pull request #7566 from wanghaox/iou_sim
add iou similarity operator
2 parents 9536c4e + fa10f03 commit d43932c

File tree

4 files changed

+262
-0
lines changed

4 files changed

+262
-0
lines changed

paddle/operators/iou_similarity_op.cc

Lines changed: 96 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,96 @@
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/iou_similarity_op.h"
16+
17+
namespace paddle {
18+
namespace operators {
19+
20+
class IOUSimilarityOp : 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->HasInput("X"),
27+
"Input(X) of IOUSimilarityOp should not be null.");
28+
PADDLE_ENFORCE(ctx->HasInput("Y"),
29+
"Input(Y) of IOUSimilarityOp should not be null.");
30+
auto x_dims = ctx->GetInputDim("X");
31+
auto y_dims = ctx->GetInputDim("Y");
32+
33+
PADDLE_ENFORCE_EQ(x_dims.size(), 2UL, "The rank of Input(X) must be 2.");
34+
PADDLE_ENFORCE_EQ(x_dims[1], 4UL, "The shape of X is [N, 4]");
35+
PADDLE_ENFORCE_EQ(y_dims.size(), 2UL, "The rank of Input(Y) must be 2.");
36+
PADDLE_ENFORCE_EQ(y_dims[1], 4UL, "The shape of Y is [M, 4]");
37+
38+
ctx->ShareLoD("X", /*->*/ "Out");
39+
ctx->SetOutputDim("Out", framework::make_ddim({x_dims[0], y_dims[0]}));
40+
}
41+
};
42+
43+
class IOUSimilarityOpMaker : public framework::OpProtoAndCheckerMaker {
44+
public:
45+
IOUSimilarityOpMaker(OpProto *proto, OpAttrChecker *op_checker)
46+
: OpProtoAndCheckerMaker(proto, op_checker) {
47+
AddInput("X",
48+
"(LoDTensor, default LoDTensor<float>) "
49+
"Box list X is a 2-D LoDTensor with shape [N, 4] holds N boxes, "
50+
"each box is represented as [xmin, ymin, xmax, ymax], "
51+
"the shape of X is [N, 4]. [xmin, ymin] is the left top "
52+
"coordinate of the box if the input is image feature map, they "
53+
"are close to the origin of the coordinate system. "
54+
"[xmax, ymax] is the right bottom coordinate of the box. "
55+
"This tensor can contain LoD information to represent a batch "
56+
"of inputs. One instance of this batch can contain different "
57+
"numbers of entities.");
58+
AddInput("Y",
59+
"(Tensor, default Tensor<float>) "
60+
"Box list Y holds M boxes, each box is represented as "
61+
"[xmin, ymin, xmax, ymax], the shape of X is [N, 4]. "
62+
"[xmin, ymin] is the left top coordinate of the box if the "
63+
"input is image feature map, and [xmax, ymax] is the right "
64+
"bottom coordinate of the box.");
65+
66+
AddOutput("Out",
67+
"(LoDTensor, the lod is same as input X) The output of "
68+
"iou_similarity op, a tensor with shape [N, M] "
69+
"representing pairwise iou scores.");
70+
71+
AddComment(R"DOC(
72+
IOU Similarity Operator.
73+
Computes intersection-over-union (IOU) between two box lists.
74+
Box list 'X' should be a LoDTensor and 'Y' is a common Tensor,
75+
boxes in 'Y' are shared by all instance of the batched inputs of X.
76+
Given two boxes A and B, the calculation of IOU is as follows:
77+
78+
$$
79+
IOU(A, B) =
80+
\frac{area(A\cap B)}{area(A)+area(B)-area(A\cap B)}
81+
$$
82+
83+
)DOC");
84+
}
85+
};
86+
} // namespace operators
87+
} // namespace paddle
88+
89+
namespace ops = paddle::operators;
90+
REGISTER_OP_WITHOUT_GRADIENT(iou_similarity, ops::IOUSimilarityOp,
91+
ops::IOUSimilarityOpMaker);
92+
93+
REGISTER_OP_CPU_KERNEL(
94+
iou_similarity,
95+
ops::IOUSimilarityKernel<paddle::platform::CPUDeviceContext, float>,
96+
ops::IOUSimilarityKernel<paddle::platform::CPUDeviceContext, double>);

paddle/operators/iou_similarity_op.cu

Lines changed: 21 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,21 @@
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/iou_similarity_op.h"
16+
17+
namespace ops = paddle::operators;
18+
REGISTER_OP_CUDA_KERNEL(
19+
iou_similarity,
20+
ops::IOUSimilarityKernel<paddle::platform::CUDADeviceContext, float>,
21+
ops::IOUSimilarityKernel<paddle::platform::CUDADeviceContext, double>);

paddle/operators/iou_similarity_op.h

Lines changed: 90 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,90 @@
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/platform/for_range.h"
18+
19+
template <typename T>
20+
inline HOSTDEVICE T IOUSimilarity(T xmin1, T ymin1, T xmax1, T ymax1, T xmin2,
21+
T ymin2, T xmax2, T ymax2) {
22+
constexpr T zero = static_cast<T>(0);
23+
T area1 = (ymax1 - ymin1) * (xmax1 - xmin1);
24+
T area2 = (ymax2 - ymin2) * (xmax2 - xmin2);
25+
T inter_xmax = xmax1 > xmax2 ? xmax2 : xmax1;
26+
T inter_ymax = ymax1 > ymax2 ? ymax2 : ymax1;
27+
T inter_xmin = xmin1 > xmin2 ? xmin1 : xmin2;
28+
T inter_ymin = ymin1 > ymin2 ? ymin1 : ymin2;
29+
T inter_height = inter_ymax - inter_ymin;
30+
T inter_width = inter_xmax - inter_xmin;
31+
inter_height = inter_height > zero ? inter_height : zero;
32+
inter_width = inter_width > zero ? inter_width : zero;
33+
T inter_area = inter_width * inter_height;
34+
T union_area = area1 + area2 - inter_area;
35+
T sim_score = inter_area / union_area;
36+
return sim_score;
37+
}
38+
39+
template <typename T>
40+
struct IOUSimilarityFunctor {
41+
IOUSimilarityFunctor(const T* x, const T* y, T* z, int cols)
42+
: x_(x), y_(y), z_(z), cols_(static_cast<size_t>(cols)) {}
43+
44+
inline HOSTDEVICE void operator()(size_t row_id) const {
45+
T x_min1 = x_[row_id * 4];
46+
T y_min1 = x_[row_id * 4 + 1];
47+
T x_max1 = x_[row_id * 4 + 2];
48+
T y_max1 = x_[row_id * 4 + 3];
49+
for (size_t i = 0; i < cols_; ++i) {
50+
T x_min2 = y_[i * 4];
51+
T y_min2 = y_[i * 4 + 1];
52+
T x_max2 = y_[i * 4 + 2];
53+
T y_max2 = y_[i * 4 + 3];
54+
55+
T sim = IOUSimilarity(x_min1, y_min1, x_max1, y_max1, x_min2, y_min2,
56+
x_max2, y_max2);
57+
58+
z_[row_id * cols_ + i] = sim;
59+
}
60+
}
61+
const T* x_;
62+
const T* y_;
63+
T* z_;
64+
const size_t cols_;
65+
};
66+
67+
namespace paddle {
68+
namespace operators {
69+
70+
template <typename DeviceContext, typename T>
71+
class IOUSimilarityKernel : public framework::OpKernel<T> {
72+
public:
73+
void Compute(const framework::ExecutionContext& ctx) const override {
74+
const framework::LoDTensor* in_x = ctx.Input<framework::LoDTensor>("X");
75+
const framework::Tensor* in_y = ctx.Input<framework::Tensor>("Y");
76+
framework::LoDTensor* out = ctx.Output<framework::LoDTensor>("Out");
77+
78+
int x_n = in_x->dims()[0];
79+
int y_n = in_y->dims()[0];
80+
IOUSimilarityFunctor<T> functor(in_x->data<T>(), in_y->data<T>(),
81+
out->mutable_data<T>(ctx.GetPlace()), y_n);
82+
83+
platform::ForRange<DeviceContext> for_range(
84+
static_cast<const DeviceContext&>(ctx.device_context()), x_n);
85+
for_range(functor);
86+
}
87+
}; // namespace operators
88+
89+
} // namespace operators
90+
} // namespace paddle
Lines changed: 55 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,55 @@
1+
# Copyright (c) 2018 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+
import unittest
16+
import numpy as np
17+
import sys
18+
import math
19+
from op_test import OpTest
20+
21+
22+
class TestIOUSimilarityOp(OpTest):
23+
def test_check_output(self):
24+
self.check_output()
25+
26+
def setUp(self):
27+
self.op_type = "iou_similarity"
28+
self.boxes1 = np.array(
29+
[[4.0, 3.0, 7.0, 5.0], [5.0, 6.0, 10.0, 7.0]]).astype('float32')
30+
self.boxes2 = np.array([[3.0, 4.0, 6.0, 8.0], [14.0, 14.0, 15.0, 15.0],
31+
[0.0, 0.0, 20.0, 20.0]]).astype('float32')
32+
self.output = np.array(
33+
[[2.0 / 16.0, 0, 6.0 / 400.0],
34+
[1.0 / 16.0, 0.0, 5.0 / 400.0]]).astype('float32')
35+
36+
self.inputs = {'X': self.boxes1, 'Y': self.boxes2}
37+
38+
self.outputs = {'Out': self.output}
39+
40+
41+
class TestIOUSimilarityOpWithLoD(TestIOUSimilarityOp):
42+
def test_check_output(self):
43+
self.check_output()
44+
45+
def setUp(self):
46+
super(TestIOUSimilarityOpWithLoD, self).setUp()
47+
self.boxes1_lod = [[0, 1, 2]]
48+
self.output_lod = [[0, 1, 2]]
49+
50+
self.inputs = {'X': (self.boxes1, self.boxes1_lod), 'Y': self.boxes2}
51+
self.outputs = {'Out': (self.output, self.output_lod)}
52+
53+
54+
if __name__ == '__main__':
55+
unittest.main()

0 commit comments

Comments
 (0)