Skip to content
Merged
Show file tree
Hide file tree
Changes from all commits
Commits
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
28 changes: 28 additions & 0 deletions paddle/phi/kernels/funcs/detection/bbox_util.h
Original file line number Diff line number Diff line change
@@ -0,0 +1,28 @@
/* Copyright (c) 2022 PaddlePaddle Authors. All Rights Reserved.

Licensed under the Apache License, Version 2.0 (the "License");
you may not use this file except in compliance with the License.
You may obtain a copy of the License at

http://www.apache.org/licenses/LICENSE-2.0

Unless required by applicable law or agreed to in writing, software
distributed under the License is distributed on an "AS IS" BASIS,
WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
See the License for the specific language governing permissions and
limitations under the License. */

#pragma once

namespace phi {
namespace funcs {

struct RangeInitFunctor {
int start_;
int delta_;
int *out_;
__device__ void operator()(size_t i) { out_[i] = start_ + i * delta_; }
};

} // namespace funcs
} // namespace phi
17 changes: 14 additions & 3 deletions paddle/phi/kernels/gpu/distribute_fpn_proposals_kernel.cu
Original file line number Diff line number Diff line change
Expand Up @@ -24,14 +24,14 @@ namespace cub = hipcub;

#include "paddle/phi/backends/gpu/gpu_context.h"
#include "paddle/phi/core/kernel_registry.h"
#include "paddle/phi/kernels/funcs/detection/bbox_util.h"
#include "paddle/phi/kernels/funcs/distribute_fpn_proposals_functor.h"
#include "paddle/phi/kernels/funcs/for_range.h"
#include "paddle/phi/kernels/funcs/gather.cu.h"
#include "paddle/phi/kernels/funcs/math_function.h"

#include "paddle/fluid/memory/allocation/allocator.h"
#include "paddle/fluid/memory/memcpy.h"
#include "paddle/fluid/operators/detection/bbox_util.h"
#include "paddle/phi/backends/gpu/gpu_primitives.h"

namespace phi {
Expand Down Expand Up @@ -62,7 +62,18 @@ __global__ void GPUDistFpnProposalsHelper(const int nthreads,
const T* offset_roi = rois + i * BBoxSize;
int roi_batch_ind = roi_batch_id_data[i];
// get the target level of current rois
T roi_area = paddle::operators::RoIArea(offset_roi, pixel_offset);
T roi_area;
if (offset_roi[2] < offset_roi[0] || offset_roi[3] < offset_roi[1]) {
roi_area = static_cast<T>(0.);
} else {
const T w = offset_roi[2] - offset_roi[0];
const T h = offset_roi[3] - offset_roi[1];
if (pixel_offset) {
roi_area = (w + 1) * (h + 1);
} else {
roi_area = w * h;
}
}
T roi_scale = sqrt(roi_area);
int tgt_lvl = floor(
log2(roi_scale / static_cast<T>(refer_scale) + (T)1e-8) + refer_level);
Expand Down Expand Up @@ -155,7 +166,7 @@ void DistributeFpnProposalsKernel(
index_in_t.Resize({roi_num});
int* idx_in = dev_ctx.template Alloc<int>(&index_in_t);
funcs::ForRange<phi::GPUContext> for_range(dev_ctx, roi_num);
for_range(paddle::operators::RangeInitFunctor{0, 1, idx_in});
for_range(funcs::RangeInitFunctor{0, 1, idx_in});

DenseTensor keys_out_t;
keys_out_t.Resize({roi_num});
Expand Down
10 changes: 2 additions & 8 deletions paddle/phi/kernels/gpu/generate_proposals_kernel.cu
Original file line number Diff line number Diff line change
Expand Up @@ -26,6 +26,7 @@ namespace cub = hipcub;

#include "paddle/phi/backends/gpu/gpu_context.h"
#include "paddle/phi/core/kernel_registry.h"
#include "paddle/phi/kernels/funcs/detection/bbox_util.h"
#include "paddle/phi/kernels/funcs/for_range.h"
#include "paddle/phi/kernels/funcs/gather.cu.h"
#include "paddle/phi/kernels/funcs/math_function.h"
Expand All @@ -38,13 +39,6 @@ int const kThreadsPerBlock = sizeof(uint64_t) * 8;

static const double kBBoxClipDefault = std::log(1000.0 / 16.0);

struct RangeInitFunctor {
int start_;
int delta_;
int *out_;
__device__ void operator()(size_t i) { out_[i] = start_ + i * delta_; }
};

template <typename T>
static void SortDescending(const phi::GPUContext &ctx,
const DenseTensor &value,
Expand All @@ -55,7 +49,7 @@ static void SortDescending(const phi::GPUContext &ctx,
index_in_t.Resize(phi::make_ddim({num}));
int *idx_in = ctx.template Alloc<int>(&index_in_t);
phi::funcs::ForRange<phi::GPUContext> for_range(ctx, num);
for_range(RangeInitFunctor{0, 1, idx_in});
for_range(funcs::RangeInitFunctor{0, 1, idx_in});

index_out->Resize(phi::make_ddim({num}));
int *idx_out = ctx.template Alloc<int>(index_out);
Expand Down