180 lines
6.9 KiB
C++
180 lines
6.9 KiB
C++
/* Copyright 2021 The TensorFlow 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.
|
|
|
|
==============================================================================*/
|
|
|
|
#include "llvm/ADT/STLExtras.h"
|
|
#include "llvm/ADT/SmallSet.h"
|
|
#include "llvm/ADT/SmallVector.h"
|
|
#include "mlir-hlo/Dialect/mhlo/IR/chlo_ops.h"
|
|
#include "mlir-hlo/Dialect/mhlo/IR/hlo_ops.h"
|
|
#include "mlir-hlo/Dialect/mhlo/transforms/passes.h"
|
|
#include "mlir-hlo/Dialect/mhlo/transforms/rewriters.h"
|
|
#include "mlir/Dialect/StandardOps/IR/Ops.h"
|
|
#include "mlir/Dialect/Tensor/IR/Tensor.h"
|
|
#include "mlir/IR/BlockAndValueMapping.h"
|
|
#include "mlir/IR/BuiltinOps.h"
|
|
#include "mlir/IR/BuiltinTypes.h"
|
|
#include "mlir/IR/Operation.h"
|
|
#include "mlir/IR/PatternMatch.h"
|
|
#include "mlir/Interfaces/InferTypeOpInterface.h"
|
|
#include "mlir/Pass/Pass.h"
|
|
#include "mlir/Transforms/GreedyPatternRewriteDriver.h"
|
|
|
|
namespace mlir {
|
|
|
|
// Needed to build `llvm::SmallSet`s of `mlir::Value`s.
|
|
static bool operator<(const Value &lhs, const Value &rhs) {
|
|
return lhs.getAsOpaquePointer() < rhs.getAsOpaquePointer();
|
|
}
|
|
|
|
namespace mhlo {
|
|
namespace {
|
|
|
|
/// Identify clusters of operations that can be rank-specialized together. The
|
|
/// required traits for clustered operations are:
|
|
/// - Element-wise: All operations in the group must be element-wise. This
|
|
/// allows to reshape operands before applying the operations as well as
|
|
/// reshaping the result to the desired shape afterwards. This way, we can,
|
|
/// e.g., apply unary ops to a completely flattened operand and restore the
|
|
/// original shape afterwards.
|
|
/// - Broadcasting semantics: All operations must implement broadcasting
|
|
/// semantics. Most importantly, this allows extending operand shapes such
|
|
/// that they match in rank.
|
|
/// - Shape reification: All operations must implement
|
|
/// `InferShapedTypeOpInterface`. This is later needed to compute and to
|
|
/// restore the desired result shape.
|
|
|
|
bool IsClusterable(Operation *op) {
|
|
if (!llvm::isa<InferShapedTypeOpInterface>(op)) return false;
|
|
unsigned int num_operands = op->getNumOperands();
|
|
if (num_operands == 0) return false;
|
|
if (num_operands == 1) return op->hasTrait<OpTrait::Elementwise>();
|
|
return op->hasTrait<chlo::OpTrait::BroadcastingElementwise>() &&
|
|
op->hasTrait<chlo::OpTrait::Broadcasting>();
|
|
}
|
|
|
|
struct RankSpecializationClusterPattern : public RewritePattern {
|
|
explicit RankSpecializationClusterPattern(MLIRContext *ctx)
|
|
: RewritePattern(MatchAnyOpTypeTag(), /*benefit=*/1, ctx) {}
|
|
|
|
LogicalResult matchAndRewrite(Operation *root_op,
|
|
PatternRewriter &rewriter) const override {
|
|
// Only apply to operations that have not been clustered yet.
|
|
if (root_op->getParentOfType<chlo::RankSpecializationClusterOp>()) {
|
|
return failure();
|
|
}
|
|
|
|
// Only cluster when rank specialization is needed.
|
|
if (!IsClusterable(root_op) ||
|
|
!llvm::any_of(root_op->getOperandTypes(),
|
|
[](Type ty) { return ty.isa<UnrankedTensorType>(); })) {
|
|
return failure();
|
|
}
|
|
|
|
// Collect all collectively rank specializable ops.
|
|
SmallVector<Operation *, 16> cluster;
|
|
llvm::SmallSet<Value, 16> operand_set;
|
|
llvm::SmallSet<Value, 16> result_set;
|
|
Operation *new_op = root_op;
|
|
while (new_op != nullptr && IsClusterable(new_op)) {
|
|
// Find results that escape the cluster.
|
|
for (OpOperand &use : new_op->getUses()) {
|
|
if (!llvm::is_contained(cluster, use.getOwner()))
|
|
result_set.insert(use.get());
|
|
}
|
|
|
|
// Update cluster operands.
|
|
for (OpResult v : new_op->getResults()) operand_set.erase(Value(v));
|
|
for (OpOperand &v : new_op->getOpOperands()) operand_set.insert(v.get());
|
|
|
|
cluster.push_back(new_op);
|
|
new_op = new_op->getPrevNode();
|
|
}
|
|
|
|
// Create `RankSpecializationClusterOp`.
|
|
auto operands = llvm::to_vector<16>(operand_set);
|
|
auto results = llvm::to_vector<16>(result_set);
|
|
auto result_types = llvm::to_vector<16>(
|
|
llvm::map_range(result_set, [](Value v) { return v.getType(); }));
|
|
Location loc = root_op->getLoc();
|
|
auto cluster_op = rewriter.create<chlo::RankSpecializationClusterOp>(
|
|
loc, result_types, operands);
|
|
|
|
// Create body block.
|
|
auto operand_types = llvm::to_vector<16>(
|
|
llvm::map_range(operand_set, [](Value v) { return v.getType(); }));
|
|
Block *block = rewriter.createBlock(&cluster_op.body(), {}, operand_types);
|
|
|
|
// Copy operations into the body.
|
|
BlockAndValueMapping bvm;
|
|
for (auto it : llvm::zip(operands, block->getArguments()))
|
|
bvm.map(std::get<0>(it), std::get<1>(it));
|
|
rewriter.setInsertionPointToStart(block);
|
|
for (Operation *it : llvm::reverse(cluster)) rewriter.clone(*it, bvm);
|
|
|
|
// Create `RankSpecializationClusterYieldOp`.
|
|
auto mapped_results = llvm::to_vector<16>(
|
|
llvm::map_range(results, [&](Value v) { return bvm.lookup(v); }));
|
|
rewriter.create<chlo::RankSpecializationClusterYieldOp>(loc,
|
|
mapped_results);
|
|
|
|
// Replace original ops with the new results.
|
|
for (auto it : llvm::zip(results, cluster_op.results()))
|
|
bvm.map(std::get<0>(it), std::get<1>(it));
|
|
for (Operation *it : cluster) {
|
|
if (it->getUses().empty()) {
|
|
rewriter.eraseOp(it);
|
|
continue;
|
|
}
|
|
auto replacements = llvm::to_vector<16>(llvm::map_range(
|
|
it->getResults(), [&](Value v) { return bvm.lookup(v); }));
|
|
rewriter.replaceOp(root_op, replacements);
|
|
}
|
|
|
|
return success();
|
|
}
|
|
};
|
|
|
|
struct RankSpecializationClusterPass
|
|
: public PassWrapper<RankSpecializationClusterPass, FunctionPass> {
|
|
void getDependentDialects(DialectRegistry ®istry) const override {
|
|
registry.insert<mhlo::MhloDialect, chlo::HloClientDialect>();
|
|
}
|
|
|
|
void runOnFunction() override {
|
|
MLIRContext *ctx = &getContext();
|
|
RewritePatternSet patterns(ctx);
|
|
mhlo::PopulateRankSpecializationClusterPatterns(ctx, &patterns);
|
|
if (failed(
|
|
applyPatternsAndFoldGreedily(getFunction(), std::move(patterns)))) {
|
|
return signalPassFailure();
|
|
}
|
|
}
|
|
};
|
|
|
|
} // namespace
|
|
|
|
void PopulateRankSpecializationClusterPatterns(
|
|
MLIRContext *context, OwningRewritePatternList *patterns) {
|
|
patterns->insert<RankSpecializationClusterPattern>(context);
|
|
}
|
|
|
|
std::unique_ptr<FunctionPass> createRankSpecializationClusterPass() {
|
|
return std::make_unique<RankSpecializationClusterPass>();
|
|
}
|
|
|
|
} // namespace mhlo
|
|
} // namespace mlir
|