Skip to content
Merged
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
18 changes: 18 additions & 0 deletions aten/src/ATen/native/cuda/Reduce.cuh
Original file line number Diff line number Diff line change
Expand Up @@ -822,6 +822,23 @@ struct ReduceOp {
} else {
index_t input_offset = threadIdx.y;
index_t step = blockDim.y;
#ifdef USE_ROCM // Prefetch loads to better hide their latency
#define PRFCH 4
for (; input_offset < config.ctas_per_output; input_offset += step*PRFCH) {
arg_vec_t next[PRFCH];
#pragma unroll
for (int u = 0; (u < PRFCH) && (input_offset + u*step < config.ctas_per_output); u++) {
index_t idx = config.staging_memory_offset(input_offset + u*step);
next[u] = reduce_buffer[idx];
}
for (int u = 0; (u < PRFCH) && (input_offset + u*step < config.ctas_per_output); u++) {
#pragma unroll
for (int i = 0; i < output_vec_size; i++) {
value[i] = ops.combine(value[i], next[u][i]);
}
}
}
#else
for (; input_offset < config.ctas_per_output; input_offset += step) {
index_t idx = config.staging_memory_offset(input_offset);
arg_vec_t next = reduce_buffer[idx];
Expand All @@ -830,6 +847,7 @@ struct ReduceOp {
value[i] = ops.combine(value[i], next[i]);
}
}
#endif
}
value = block_y_reduce<output_vec_size>(value, shared_memory);
if (config.should_block_x_reduce()) {
Expand Down