lrn_op.cu 6.1 KB
Newer Older
G
gongweibao 已提交
1 2 3 4 5 6 7 8 9 10 11 12 13 14 15 16
/* Copyright (c) 2016 PaddlePaddle Authors. All Rights Reserve.

   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. */

#include "paddle/operators/lrn_op.h"

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
namespace paddle {
namespace operators {

template <typename T>
__global__ void KeCMRNormFillScale(int img_size, const T* in, T* mid, int C,
                                   int H, int W, int size, T k, T alpha) {
  const int idx = threadIdx.x + blockIdx.x * blockDim.x;
  if (idx < img_size) {
    const int w = idx % W;
    const int h = (idx / W) % H;
    const int n = idx / W / H;
    const int offset = (n * C * H + h) * W + w;

    in += offset;
    mid += offset;
    const int step = H * W;
    const int pre_pad = (size - 1) / 2;
    const int post_pad = size - pre_pad - 1;

    T accum = 0;
    int index = 0;
    while (index < C + post_pad) {
      if (index < C) {
        T val = in[index * step];
        accum += val * val;
      }
      if (index >= size) {
        T val = in[(index - size) * step];
        accum -= val * val;
      }
      if (index >= post_pad) {
        mid[(index - post_pad) * step] = k + accum * alpha;
      }
      ++index;
    }
  }
}

template <typename T>
__global__ void KeCMRNormOutput(int input_size, const T* in, const T* mid,
                                T negative_beta, T* out) {
  const int index = threadIdx.x + blockIdx.x * blockDim.x;
  if (index < input_size) {
    out[index] = in[index] * pow(mid[index], negative_beta);
  }
}

template <typename T>
void CrossMapNormal(const framework::ExecutionContext& ctx, const T* inputs,
                    T* outputs, T* mid, int N, int C, int H, int W, int n, T k,
                    T alpha, T beta) {
  int img_size = N * H * W;
  const int block_size = 1024;
  int grid_size = (img_size + block_size - 1) / block_size;

  KeCMRNormFillScale<
      T><<<grid_size, block_size, 0, ctx.cuda_device_context().stream()>>>(
      img_size, inputs, mid, C, H, W, n, k, alpha);

  int input_size = N * H * W * C;
  grid_size = (input_size + block_size - 1) / block_size;
  KeCMRNormOutput<
      T><<<grid_size, block_size, 0, ctx.cuda_device_context().stream()>>>(
      input_size, inputs, mid, -beta, outputs);
}

template <typename T>
struct LRNFunctor<platform::GPUPlace, T> {
  void operator()(const framework::ExecutionContext& ctx,
                  const framework::Tensor& input, framework::Tensor* out,
                  framework::Tensor* mid, int N, int C, int H, int W, int n,
                  T k, T alpha, T beta) {
    CrossMapNormal<T>(
        ctx, input.data<T>(), out->mutable_data<T>(ctx.GetPlace()),
        mid->mutable_data<T>(ctx.GetPlace()), N, C, H, W, n, k, alpha, beta);
  }
};

template struct LRNFunctor<platform::GPUPlace, float>;
template struct LRNFunctor<platform::GPUPlace, double>;
G
gongweibao 已提交
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
template <typename T>
__global__ void KeCMRNormDiff(int img_size, const T* x, const T* out,
                              const T* mid, T* x_g, const T* out_g, int C,
                              int H, int W, int size, T negative_beta,
                              T ratio) {
  const int idx = threadIdx.x + blockIdx.x * blockDim.x;
  if (idx < img_size) {
    const int w = idx % W;
    const int h = (idx / W) % H;
    const int n = idx / W / H;
    const int offset = (n * C * H + h) * W + w;
    x += offset;
    out += offset;
    mid += offset;
    out_g += offset;
    x_g += offset;

    const int step = H * W;
    const int pre_pad = size - (size + 1) / 2;
    const int post_pad = size - pre_pad - 1;

    int index = 0;
    T accum = 0;
    // TODO(gongwb): optimize this with thread shared array.
    while (index < C + post_pad) {
      if (index < C) {
        x_g[index * step] = 0.0;
        accum += out_g[index * step] * out[index * step] / mid[index * step];
      }
      if (index >= size) {
        accum -= out_g[(index - size) * step] * out[(index - size) * step] /
                 mid[(index - size) * step];
      }
      if (index >= post_pad) {
        x_g[(index - post_pad) * step] +=
            out_g[(index - post_pad) * step] *
                pow(mid[(index - post_pad) * step], negative_beta) -
            ratio * x[(index - post_pad) * step] * accum;
      }
      ++index;
    }
  }
}

template <typename T>
void CrossMapNormalGrad(const framework::ExecutionContext& ctx, const T* x,
                        const T* out, const T* mid, T* x_g, const T* out_g,
                        int N, int C, int H, int W, int n, T alpha, T beta) {
  int img_size = N * H * W;

  const int block_size = 1024;
  int grid_size = (img_size + block_size - 1) / block_size;

  KeCMRNormDiff<
      T><<<grid_size, block_size, 0, ctx.cuda_device_context().stream()>>>(
      img_size, x, out, mid, x_g, out_g, C, H, W, n, -beta,
      2.0f * alpha * beta);
}

template <typename T>
struct LRNGradFunctor<platform::GPUPlace, T> {
  void operator()(const framework::ExecutionContext& ctx,
                  const framework::Tensor& x, const framework::Tensor& out,
                  const framework::Tensor& mid, framework::Tensor* x_g,
                  const framework::Tensor& out_g, int N, int C, int H, int W,
                  int n, T alpha, T beta) {
    CrossMapNormalGrad<T>(ctx, x.data<T>(), out.data<T>(), mid.data<T>(),
                          x_g->mutable_data<T>(ctx.GetPlace()), out_g.data<T>(),
                          N, C, H, W, n, alpha, beta);
  }
};

template struct LRNGradFunctor<platform::GPUPlace, float>;
template struct LRNGradFunctor<platform::GPUPlace, double>;
}  // namespace operators
}  // namespace paddle

namespace ops = paddle::operators;
G
gongweibao 已提交
176 177 178
REGISTER_OP_GPU_KERNEL(lrn, ops::LRNKernel<paddle::platform::GPUPlace, float>);
REGISTER_OP_GPU_KERNEL(lrn_grad,
                       ops::LRNGradKernel<paddle::platform::GPUPlace, float>);