malloc_test.cu 7.0 KB
Newer Older
1 2 3 4 5 6 7 8 9 10 11 12 13 14
// Copyright (c) 2018 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.

15 16 17 18 19 20 21 22 23
#include <thread>  // NOLINT
#include <vector>

#include "gtest/gtest.h"
#include "paddle/fluid/memory/allocation/allocator_facade.h"
#include "paddle/fluid/memory/malloc.h"
#include "paddle/fluid/platform/device_context.h"
#include "paddle/phi/core/stream.h"

24
#ifdef PADDLE_WITH_CUDA
25 26
#include <cuda.h>
#include <cuda_runtime.h>
27 28 29 30 31 32
#endif

#ifdef PADDLE_WITH_HIP
#include <hip/hip_runtime.h>
#endif

33 34 35 36 37 38 39 40 41 42 43 44 45 46 47 48 49 50 51
namespace paddle {
namespace memory {

const int NUM_STREAMS = 8;
const int N = 2;
const float DELTA = 1e-1;

using CudaDevCtxVec = std::vector<std::unique_ptr<platform::CUDADeviceContext>>;

__global__ void kernel(float *x, int n) {
  int tid = threadIdx.x + blockIdx.x * blockDim.x;
  for (int i = tid; i < n; i += blockDim.x * gridDim.x) {
    x[i] = 3.14159 * i;
  }
}

void CheckKernelOutput(float *x, int n) {
  auto host_x = std::unique_ptr<float[]>(new float[n]);
  for (int i = 0; i < n; ++i) {
52
#ifdef PADDLE_WITH_HIP
53 54 55
    EXPECT_TRUE(
        hipSuccess ==
        hipMemcpy(host_x.get(), x, n * sizeof(float), hipMemcpyDeviceToHost));
56
#else
57 58 59
    EXPECT_TRUE(
        cudaSuccess ==
        cudaMemcpy(host_x.get(), x, n * sizeof(float), cudaMemcpyDeviceToHost));
60
#endif
61 62 63 64 65
    EXPECT_GE(host_x[i] + DELTA, 3.14159f * i);
    EXPECT_LE(host_x[i] - DELTA, 3.14159f * i);
  }
}

66 67
void MultiStreamCompute(float **data,
                        float **second_data,
68 69 70 71 72
                        const platform::CUDADeviceContext &ctx) {
  // multi-streams
  AllocationPtr allocation_ptr = Alloc(ctx, N * sizeof(float));
  EXPECT_GE(allocation_ptr->size(), N * sizeof(float));
  *data = reinterpret_cast<float *>(allocation_ptr->ptr());
73 74 75
#ifdef PADDLE_WITH_HIP
  hipLaunchKernelGGL((kernel), dim3(1), dim3(64), 0, ctx.stream(), *data, N);
#else
76
  kernel<<<1, 64, 0, ctx.stream()>>>(*data, N);
77
#endif
78 79 80 81 82

  // allocate and compute on same stream again
  allocation_ptr = Alloc(ctx, N * sizeof(float));
  EXPECT_GE(allocation_ptr->size(), N * sizeof(float));
  *second_data = reinterpret_cast<float *>(allocation_ptr->ptr());
83
#ifdef PADDLE_WITH_HIP
84 85
  hipLaunchKernelGGL(
      (kernel), dim3(1), dim3(64), 0, ctx.stream(), *second_data, N);
86
#else
87
  kernel<<<1, 64, 0, ctx.stream()>>>(*second_data, N);
88
#endif
89 90 91 92
}

TEST(Malloc, CUDADeviceContextMultiStream) {
  auto place = platform::CUDAPlace(0);
L
Leo Chen 已提交
93
  platform::SetDeviceId(0);
94 95 96 97 98 99 100 101 102 103

  AllocationPtr main_stream_alloc_ptr = Alloc(place, N * sizeof(float));
  EXPECT_GE(main_stream_alloc_ptr->size(), N * sizeof(float));
  float *main_stream_data =
      reinterpret_cast<float *>(main_stream_alloc_ptr->ptr());

  float *data[NUM_STREAMS];
  float *second_data[NUM_STREAMS];
  CudaDevCtxVec dev_ctx;

104 105 106 107
// default stream
#ifdef PADDLE_WITH_HIP
  hipLaunchKernelGGL((kernel), dim3(1), dim3(64), 0, 0, main_stream_data, N);
#else
108
  kernel<<<1, 64>>>(main_stream_data, N);
109
#endif
110 111 112
  main_stream_alloc_ptr.reset();

  for (int i = 0; i < NUM_STREAMS; ++i) {
W
Wilber 已提交
113 114 115 116 117 118 119 120 121 122 123 124 125
    auto ctx = std::unique_ptr<platform::CUDADeviceContext>(
        new platform::CUDADeviceContext(place));
    ctx->SetAllocator(paddle::memory::allocation::AllocatorFacade::Instance()
                          .GetAllocator(place, ctx->stream())
                          .get());
    ctx->SetHostAllocator(
        paddle::memory::allocation::AllocatorFacade::Instance()
            .GetAllocator(paddle::platform::CPUPlace())
            .get());
    ctx->SetZeroAllocator(
        paddle::memory::allocation::AllocatorFacade::Instance()
            .GetZeroAllocator(place)
            .get());
W
wanghuancoder 已提交
126 127 128 129
    ctx->SetPinnedAllocator(
        paddle::memory::allocation::AllocatorFacade::Instance()
            .GetAllocator(paddle::platform::CUDAPinnedPlace())
            .get());
W
Wilber 已提交
130 131
    ctx->PartialInitWithAllocator();
    dev_ctx.emplace_back(std::move(ctx));
132 133 134
    MultiStreamCompute(&data[i], &second_data[i], *dev_ctx[i]);
  }

135 136 137
#ifdef PADDLE_WITH_HIP
  EXPECT_TRUE(hipSuccess == hipDeviceSynchronize());
#else
138
  EXPECT_TRUE(cudaSuccess == cudaDeviceSynchronize());
139
#endif
140 141 142 143 144 145 146 147
  for (int i = 0; i < NUM_STREAMS; ++i) {
    CheckKernelOutput(data[i], N);
    CheckKernelOutput(second_data[i], N);
  }
}

TEST(Malloc, CUDADeviceContextMultiThreadMultiStream) {
  auto place = platform::CUDAPlace(0);
L
Leo Chen 已提交
148
  platform::SetDeviceId(0);
149 150 151 152 153 154 155 156 157 158 159

  AllocationPtr main_stream_alloc_ptr = Alloc(place, N * sizeof(float));
  EXPECT_GE(main_stream_alloc_ptr->size(), N * sizeof(float));
  float *main_stream_data =
      reinterpret_cast<float *>(main_stream_alloc_ptr->ptr());

  float *data[NUM_STREAMS];
  float *second_data[NUM_STREAMS];
  CudaDevCtxVec dev_ctx;
  std::vector<std::thread> threads;

160 161 162 163
// default stream
#ifdef PADDLE_WITH_HIP
  hipLaunchKernelGGL((kernel), dim3(1), dim3(64), 0, 0, main_stream_data, N);
#else
164
  kernel<<<1, 64>>>(main_stream_data, N);
165
#endif
166 167 168
  main_stream_alloc_ptr.reset();

  for (int i = 0; i < NUM_STREAMS; ++i) {
W
Wilber 已提交
169 170 171 172 173 174 175 176 177 178 179 180 181
    auto ctx = std::unique_ptr<platform::CUDADeviceContext>(
        new platform::CUDADeviceContext(place));
    ctx->SetAllocator(paddle::memory::allocation::AllocatorFacade::Instance()
                          .GetAllocator(place, ctx->stream())
                          .get());
    ctx->SetHostAllocator(
        paddle::memory::allocation::AllocatorFacade::Instance()
            .GetAllocator(paddle::platform::CPUPlace())
            .get());
    ctx->SetZeroAllocator(
        paddle::memory::allocation::AllocatorFacade::Instance()
            .GetZeroAllocator(place)
            .get());
W
wanghuancoder 已提交
182 183 184 185
    ctx->SetPinnedAllocator(
        paddle::memory::allocation::AllocatorFacade::Instance()
            .GetAllocator(paddle::platform::CUDAPinnedPlace())
            .get());
W
Wilber 已提交
186 187
    ctx->PartialInitWithAllocator();
    dev_ctx.emplace_back(std::move(ctx));
188 189
    threads.push_back(std::thread(
        MultiStreamCompute, &data[i], &second_data[i], std::cref(*dev_ctx[i])));
190 191 192 193 194
  }

  for (int i = 0; i < NUM_STREAMS; ++i) {
    threads[i].join();
  }
195 196 197
#ifdef PADDLE_WITH_HIP
  EXPECT_TRUE(hipSuccess == hipDeviceSynchronize());
#else
198
  EXPECT_TRUE(cudaSuccess == cudaDeviceSynchronize());
199
#endif
200 201 202 203 204 205 206 207 208 209 210
  for (int i = 0; i < NUM_STREAMS; ++i) {
    CheckKernelOutput(data[i], N);
    CheckKernelOutput(second_data[i], N);
  }
}

TEST(Malloc, AllocZero) {
  auto place = platform::CUDAPlace(0);
  AllocationPtr allocation_ptr = Alloc(place, 0);
  EXPECT_GE(allocation_ptr->size(), 0);
}
211 212 213 214 215 216 217

TEST(Malloc, AllocWithStream) {
  size_t size = 1024;
  AllocationPtr allocation = Alloc(platform::CUDAPlace(), size, phi::Stream(0));
  EXPECT_EQ(allocation->size(), 1024);
}

218 219
}  // namespace memory
}  // namespace paddle